From 4904aa2aaa44a661ee58b505b90ba49c1f3d1427 Mon Sep 17 00:00:00 2001 From: Micalling <2687920886@qq.com> Date: Tue, 17 Sep 2024 13:26:46 +0800 Subject: [PATCH 1/4] =?UTF-8?q?=E3=80=90Hackathon=207th=20No.22=E3=80=91NO?= =?UTF-8?q?.22=20=E5=9C=A8=20paddle.audio.functional.get=5Fwindow=20?= =?UTF-8?q?=E4=B8=AD=E6=94=AF=E6=8C=81=20bartlett=20=E3=80=81=20kaiser=20?= =?UTF-8?q?=E5=92=8C=20nuttall=20=E7=AA=97=E5=87=BD=E6=95=B0?= MIME-Version: 1.0 Content-Type: text/plain; charset=UTF-8 Content-Transfer-Encoding: 8bit --- python/paddle/audio/features/layers.py | 2 + python/paddle/audio/functional/window.py | 63 +++++++++++- test/legacy_test/test_audio_functions.py | 19 +++- test/legacy_test/test_get_window.py | 116 +++++++++++++++++++++++ 4 files changed, 193 insertions(+), 7 deletions(-) create mode 100644 test/legacy_test/test_get_window.py diff --git a/python/paddle/audio/features/layers.py b/python/paddle/audio/features/layers.py index 5f72d27d854d5..36f814865bf32 100644 --- a/python/paddle/audio/features/layers.py +++ b/python/paddle/audio/features/layers.py @@ -31,6 +31,8 @@ 'hamming', 'hann', 'kaiser', + 'bartlett', + 'nuttall', 'gaussian', 'exponential', 'triang', diff --git a/python/paddle/audio/functional/window.py b/python/paddle/audio/functional/window.py index 5ea552fc5e4c6..b75b310c9ae2f 100644 --- a/python/paddle/audio/functional/window.py +++ b/python/paddle/audio/functional/window.py @@ -55,6 +55,61 @@ def _cat(x: list[Tensor], data_type: str) -> Tensor: return paddle.concat(l) +@window_function_register.register() +def _bartlett(M: int, sym: bool = True, dtype: str = 'float64') -> Tensor: + """ + Computes the Bartlett window. + This function is consistent with scipy.signal.windows.bartlett(). + """ + if _len_guards(M): + return paddle.ones((M,), dtype=dtype) + M, needs_trunc = _extend(M, sym) + + n = paddle.arange(0, M, dtype=dtype) + M = paddle.to_tensor(M, dtype=dtype) + w = paddle.where( + paddle.less_equal(n, (M - 1) / 2.0), + 2.0 * n / (M - 1), + 2.0 - 2.0 * n / (M - 1), + ) + + return _truncate(w, needs_trunc) + + +@window_function_register.register() +def _kaiser( + M: int, beta: float, sym: bool = True, dtype: str = 'float64' +) -> Tensor: + """Compute the Kaiser window. + This function is consistent with scipy.signal.windows.kaiser(). + """ + if _len_guards(M): + return paddle.ones((M,), dtype=dtype) + M, needs_trunc = _extend(M, sym) + + beta = paddle.to_tensor(beta, dtype=dtype) + + n = paddle.arange(0, M, dtype=dtype) + M = paddle.to_tensor(M, dtype=dtype) + alpha = (M - 1) / 2.0 + w = paddle.i0( + beta * paddle.sqrt(1 - ((n - alpha) / alpha) ** 2.0) + ) / paddle.i0(beta) + + return _truncate(w, needs_trunc) + + +@window_function_register.register() +def _nuttall(M: int, sym: bool = True, dtype: str = 'float64') -> Tensor: + """Nuttall window. + This function is consistent with scipy.signal.windows.nuttall(). + """ + a = paddle.to_tensor( + [0.3635819, 0.4891775, 0.1365995, 0.0106411], dtype=dtype + ) + return _general_cosine(M, a=a, sym=sym, dtype=dtype) + + @window_function_register.register() def _acosh(x: Tensor | float) -> Tensor: if isinstance(x, float): @@ -347,7 +402,7 @@ def get_window( """Return a window of a given length and type. Args: - window (Union[str, Tuple[str, float]]): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'gaussian', 'general_gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor'. + window (Union[str, Tuple[str, float]]): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'gaussian', 'general_gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor', 'bartlett', 'kaiser', 'nuttall'. win_length (int): Number of samples. fftbins (bool, optional): If True, create a "periodic" window. Otherwise, create a "symmetric" window, for use in filter design. Defaults to True. dtype (str, optional): The data type of the return window. Defaults to 'float64'. @@ -364,17 +419,16 @@ def get_window( >>> cosine_window = paddle.audio.functional.get_window('cosine', n_fft) >>> std = 7 - >>> gaussian_window = paddle.audio.functional.get_window(('gaussian',std), n_fft) + >>> gaussian_window = paddle.audio.functional.get_window(('gaussian', std), n_fft) """ sym = not fftbins - args = () if isinstance(window, tuple): winstr = window[0] if len(window) > 1: args = window[1:] elif isinstance(window, str): - if window in ['gaussian', 'exponential']: + if window in ['gaussian', 'exponential', 'kaiser']: raise ValueError( "The '" + window + "' window needs one or " "more parameters -- pass a tuple." @@ -388,7 +442,6 @@ def get_window( winfunc = window_function_register.get('_' + winstr) except KeyError as e: raise ValueError("Unknown window type.") from e - params = (win_length, *args) kwargs = {'sym': sym} return winfunc(*params, dtype=dtype, **kwargs) diff --git a/test/legacy_test/test_audio_functions.py b/test/legacy_test/test_audio_functions.py index 9e3e42086f1cf..f1f88b00f551f 100644 --- a/test/legacy_test/test_audio_functions.py +++ b/test/legacy_test/test_audio_functions.py @@ -259,6 +259,7 @@ def test_gaussian_window_and_exception(self, n_fft: int): np.testing.assert_array_almost_equal( window_scipy_exp, window_paddle_exp.numpy(), decimal=5 ) + try: window_paddle = paddle.audio.functional.get_window("hann", -1) except ValueError: @@ -292,7 +293,14 @@ def dct(n_filters, n_input): np.testing.assert_array_almost_equal(librosa_dct, paddle_dct, decimal=5) @parameterize( - [128, 256, 512], ["hamming", "hann", "triang", "bohman"], [True, False] + [128, 256, 512], + [ + "hamming", + "hann", + "triang", + "bohman", + ], + [True, False], ) def test_stft_and_spect( self, n_fft: int, window_str: str, center_flag: bool @@ -347,7 +355,14 @@ def test_stft_and_spect( ) @parameterize( - [128, 256, 512], [64, 82], ["hamming", "hann", "triang", "bohman"] + [128, 256, 512], + [64, 82], + [ + "hamming", + "hann", + "triang", + "bohman", + ], ) def test_istft(self, n_fft: int, hop_length: int, window_str: str): if len(self.waveform.shape) == 2: # (C, T) diff --git a/test/legacy_test/test_get_window.py b/test/legacy_test/test_get_window.py new file mode 100644 index 0000000000000..4d85477d67563 --- /dev/null +++ b/test/legacy_test/test_get_window.py @@ -0,0 +1,116 @@ +# Copyright (c) 2024 PaddlePaddle Authors. All Rights Reserved. +# +# Licensed under the Apache License, Version 2.0 (the "License"); +# you may not use this file except in compliance with the License. +# You may obtain a copy of the License at +# +# http://www.apache.org/licenses/LICENSE-2.0 +# +# Unless required by applicable law or agreed to in writing, software +# distributed under the License is distributed on an "AS IS" BASIS, +# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied. +# See the License for the specific language governing permissions and +# limitations under the License. +import unittest + +import numpy as np +from scipy import signal + +import paddle +import paddle.audio +from paddle.base import core + + +class TestAudioFuncitons(unittest.TestCase): + def test_bartlett_nuttall_kaiser_window(self): + paddle.disable_static( + paddle.CUDAPlace(0) + if core.is_compiled_with_cuda() + else paddle.CPUPlace() + ) + n_fft = 1 + + window_scipy_bartlett = signal.windows.bartlett(n_fft) + window_paddle_bartlett = paddle.audio.functional.get_window( + 'bartlett', n_fft + ) + window_scipy_bartlett = paddle.to_tensor( + window_scipy_bartlett, dtype=window_paddle_bartlett.dtype + ) + paddle.allclose( + window_scipy_bartlett, + window_paddle_bartlett, + atol=0.0001, + rtol=0.0001, + ) + + window_scipy_nuttall = signal.windows.nuttall(n_fft) + window_paddle_nuttall = paddle.audio.functional.get_window( + 'nuttall', n_fft + ) + window_scipy_nuttall = paddle.to_tensor( + window_scipy_nuttall, dtype=window_paddle_nuttall.dtype + ) + paddle.allclose( + window_scipy_nuttall, + window_paddle_nuttall, + atol=0.0001, + rtol=0.0001, + ) + + window_scipy_kaiser = signal.windows.kaiser(n_fft, beta=14.0) + window_paddle_kaiser = paddle.audio.functional.get_window( + ('kaiser', 14.0), n_fft + ) + window_scipy_kaiser = paddle.to_tensor( + window_scipy_kaiser, dtype=window_paddle_kaiser.dtype + ) + paddle.allclose( + window_scipy_kaiser, window_paddle_kaiser, atol=0.0001, rtol=0.0001 + ) + + n_fft = 512 + + window_scipy_bartlett = signal.windows.bartlett(n_fft) + window_paddle_bartlett = paddle.audio.functional.get_window( + 'bartlett', n_fft + ) + window_scipy_bartlett = paddle.to_tensor( + window_scipy_bartlett, dtype=window_paddle_bartlett.dtype + ) + paddle.allclose( + window_scipy_bartlett, + window_paddle_bartlett, + atol=0.0001, + rtol=0.0001, + ) + + window_scipy_nuttall = signal.windows.nuttall(n_fft) + window_paddle_nuttall = paddle.audio.functional.get_window( + 'nuttall', n_fft + ) + window_scipy_nuttall = paddle.to_tensor( + window_scipy_nuttall, dtype=window_paddle_nuttall.dtype + ) + paddle.allclose( + window_scipy_nuttall, + window_paddle_nuttall, + atol=0.0001, + rtol=0.0001, + ) + + window_scipy_kaiser = signal.windows.kaiser(n_fft, beta=14.0) + window_paddle_kaiser = paddle.audio.functional.get_window( + ('kaiser', 14.0), n_fft + ) + window_scipy_kaiser = paddle.to_tensor( + window_scipy_kaiser, dtype=window_paddle_kaiser.dtype + ) + paddle.allclose( + window_scipy_kaiser, window_paddle_kaiser, atol=0.0001, rtol=0.0001 + ) + paddle.enable_static() + + +if __name__ == '__main__': + unittest.main() From 761b5cfdc3ac662b7717342520c2af4acfc4bc3a Mon Sep 17 00:00:00 2001 From: Micalling <2687920886@qq.com> Date: Wed, 25 Sep 2024 22:42:08 +0800 Subject: [PATCH 2/4] update --- test/legacy_test/test_get_window.py | 1 - 1 file changed, 1 deletion(-) diff --git a/test/legacy_test/test_get_window.py b/test/legacy_test/test_get_window.py index 4d85477d67563..66db869fd27aa 100644 --- a/test/legacy_test/test_get_window.py +++ b/test/legacy_test/test_get_window.py @@ -13,7 +13,6 @@ # limitations under the License. import unittest -import numpy as np from scipy import signal import paddle From 182e5ef88740fa5b787921336e0aac1606d3c9f1 Mon Sep 17 00:00:00 2001 From: Micalling <2687920886@qq.com> Date: Wed, 9 Oct 2024 10:15:44 +0800 Subject: [PATCH 3/4] update --- python/paddle/audio/features/layers.py | 8 +- test/legacy_test/test_get_window.py | 101 ++++++++++++------------- 2 files changed, 54 insertions(+), 55 deletions(-) diff --git a/python/paddle/audio/features/layers.py b/python/paddle/audio/features/layers.py index 36f814865bf32..1f578f072b8e5 100644 --- a/python/paddle/audio/features/layers.py +++ b/python/paddle/audio/features/layers.py @@ -52,7 +52,7 @@ class Spectrogram(nn.Layer): n_fft (int, optional): The number of frequency components of the discrete Fourier transform. Defaults to 512. hop_length (Optional[int], optional): The hop length of the short time FFT. If `None`, it is set to `win_length//4`. Defaults to None. win_length (Optional[int], optional): The window length of the short time FFT. If `None`, it is set to same as `n_fft`. Defaults to None. - window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'kaiser', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor'. Defaults to 'hann'. + window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor', 'bartlett', 'kaiser', 'nuttall'. Defaults to 'hann'. power (float, optional): Exponent for the magnitude spectrogram. Defaults to 2.0. center (bool, optional): Whether to pad `x` to make that the :math:`t \times hop\\_length` at the center of `t`-th frame. Defaults to True. pad_mode (str, optional): Choose padding pattern when `center` is `True`. Defaults to 'reflect'. @@ -137,7 +137,7 @@ class MelSpectrogram(nn.Layer): n_fft (int, optional): The number of frequency components of the discrete Fourier transform. Defaults to 512. hop_length (Optional[int], optional): The hop length of the short time FFT. If `None`, it is set to `win_length//4`. Defaults to None. win_length (Optional[int], optional): The window length of the short time FFT. If `None`, it is set to same as `n_fft`. Defaults to None. - window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'kaiser', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor'. Defaults to 'hann'. + window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor', 'bartlett', 'kaiser', 'nuttall'. Defaults to 'hann'. power (float, optional): Exponent for the magnitude spectrogram. Defaults to 2.0. center (bool, optional): Whether to pad `x` to make that the :math:`t \times hop\\_length` at the center of `t`-th frame. Defaults to True. pad_mode (str, optional): Choose padding pattern when `center` is `True`. Defaults to 'reflect'. @@ -244,7 +244,7 @@ class LogMelSpectrogram(nn.Layer): n_fft (int, optional): The number of frequency components of the discrete Fourier transform. Defaults to 512. hop_length (Optional[int], optional): The hop length of the short time FFT. If `None`, it is set to `win_length//4`. Defaults to None. win_length (Optional[int], optional): The window length of the short time FFT. If `None`, it is set to same as `n_fft`. Defaults to None. - window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'kaiser', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor'. Defaults to 'hann'. + window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor', 'bartlett', 'kaiser', 'nuttall'. Defaults to 'hann'. power (float, optional): Exponent for the magnitude spectrogram. Defaults to 2.0. center (bool, optional): Whether to pad `x` to make that the :math:`t \times hop\\_length` at the center of `t`-th frame. Defaults to True. pad_mode (str, optional): Choose padding pattern when `center` is `True`. Defaults to 'reflect'. @@ -352,7 +352,7 @@ class MFCC(nn.Layer): n_fft (int, optional): The number of frequency components of the discrete Fourier transform. Defaults to 512. hop_length (Optional[int], optional): The hop length of the short time FFT. If `None`, it is set to `win_length//4`. Defaults to None. win_length (Optional[int], optional): The window length of the short time FFT. If `None`, it is set to same as `n_fft`. Defaults to None. - window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'kaiser', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor'. Defaults to 'hann'. + window (str, optional): The window function applied to the signal before the Fourier transform. Supported window functions: 'hamming', 'hann', 'gaussian', 'exponential', 'triang', 'bohman', 'blackman', 'cosine', 'tukey', 'taylor', 'bartlett', 'kaiser', 'nuttall'. Defaults to 'hann'. power (float, optional): Exponent for the magnitude spectrogram. Defaults to 2.0. center (bool, optional): Whether to pad `x` to make that the :math:`t \times hop\\_length` at the center of `t`-th frame. Defaults to True. pad_mode (str, optional): Choose padding pattern when `center` is `True`. Defaults to 'reflect'. diff --git a/test/legacy_test/test_get_window.py b/test/legacy_test/test_get_window.py index 66db869fd27aa..c793ade880821 100644 --- a/test/legacy_test/test_get_window.py +++ b/test/legacy_test/test_get_window.py @@ -11,8 +11,10 @@ # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied. # See the License for the specific language governing permissions and # limitations under the License. +import itertools import unittest +from parameterized import parameterized from scipy import signal import paddle @@ -20,82 +22,80 @@ from paddle.base import core +def parameterize(*params): + return parameterized.expand(list(itertools.product(*params))) + + class TestAudioFuncitons(unittest.TestCase): - def test_bartlett_nuttall_kaiser_window(self): + def setUp(self): paddle.disable_static( paddle.CUDAPlace(0) if core.is_compiled_with_cuda() else paddle.CPUPlace() ) - n_fft = 1 - window_scipy_bartlett = signal.windows.bartlett(n_fft) - window_paddle_bartlett = paddle.audio.functional.get_window( - 'bartlett', n_fft - ) - window_scipy_bartlett = paddle.to_tensor( - window_scipy_bartlett, dtype=window_paddle_bartlett.dtype - ) + @parameterize( + [ + "hamming", + "hann", + "triang", + "bohman", + "blackman", + "cosine", + "tukey", + "taylor", + "bartlett", + "nuttall", + ], + [1, 512], + ) + def test_window(self, window_type: str, n_fft: int): + window_scipy = signal.get_window(window_type, n_fft) + window_paddle = paddle.audio.functional.get_window(window_type, n_fft) + window_scipy = paddle.to_tensor(window_scipy, dtype=window_paddle.dtype) paddle.allclose( - window_scipy_bartlett, - window_paddle_bartlett, + window_scipy, + window_paddle, atol=0.0001, rtol=0.0001, ) - window_scipy_nuttall = signal.windows.nuttall(n_fft) - window_paddle_nuttall = paddle.audio.functional.get_window( - 'nuttall', n_fft + + @parameterize([1, 512]) + def test_window_and_exception(self, n_fft: int): + window_scipy_gaussain = signal.windows.gaussian(n_fft, std=7) + window_paddle_gaussian = paddle.audio.functional.get_window( + ('gaussian', 7), n_fft, False ) - window_scipy_nuttall = paddle.to_tensor( - window_scipy_nuttall, dtype=window_paddle_nuttall.dtype + window_scipy_gaussain = paddle.to_tensor( + window_scipy_gaussain, dtype=window_paddle_gaussian.dtype ) paddle.allclose( - window_scipy_nuttall, - window_paddle_nuttall, - atol=0.0001, - rtol=0.0001, + window_scipy_gaussain, window_paddle_gaussian, atol=0.0001, rtol=0.0001 ) - window_scipy_kaiser = signal.windows.kaiser(n_fft, beta=14.0) - window_paddle_kaiser = paddle.audio.functional.get_window( - ('kaiser', 14.0), n_fft + window_scipy_general_gaussain = signal.windows.general_gaussian( + n_fft, 1, 7 ) - window_scipy_kaiser = paddle.to_tensor( - window_scipy_kaiser, dtype=window_paddle_kaiser.dtype + window_paddle_general_gaussian = paddle.audio.functional.get_window( + ('general_gaussian', 1, 7), n_fft, False ) - paddle.allclose( - window_scipy_kaiser, window_paddle_kaiser, atol=0.0001, rtol=0.0001 - ) - - n_fft = 512 - - window_scipy_bartlett = signal.windows.bartlett(n_fft) - window_paddle_bartlett = paddle.audio.functional.get_window( - 'bartlett', n_fft - ) - window_scipy_bartlett = paddle.to_tensor( - window_scipy_bartlett, dtype=window_paddle_bartlett.dtype + window_scipy_general_gaussain = paddle.to_tensor( + window_scipy_general_gaussain, dtype=window_paddle_general_gaussian.dtype ) paddle.allclose( - window_scipy_bartlett, - window_paddle_bartlett, - atol=0.0001, - rtol=0.0001, + window_scipy_gaussain, window_paddle_gaussian, atol=0.0001, rtol=0.0001 ) - window_scipy_nuttall = signal.windows.nuttall(n_fft) - window_paddle_nuttall = paddle.audio.functional.get_window( - 'nuttall', n_fft + window_scipy_exp = signal.windows.exponential(n_fft) + window_paddle_exp = paddle.audio.functional.get_window( + ('exponential', None, 1), n_fft, False ) - window_scipy_nuttall = paddle.to_tensor( - window_scipy_nuttall, dtype=window_paddle_nuttall.dtype + window_scipy_exp = paddle.to_tensor( + window_scipy_exp, dtype=window_paddle_exp.dtype ) paddle.allclose( - window_scipy_nuttall, - window_paddle_nuttall, - atol=0.0001, - rtol=0.0001, + window_scipy_exp, window_paddle_exp, atol=0.0001, rtol=0.0001 ) window_scipy_kaiser = signal.windows.kaiser(n_fft, beta=14.0) @@ -108,7 +108,6 @@ def test_bartlett_nuttall_kaiser_window(self): paddle.allclose( window_scipy_kaiser, window_paddle_kaiser, atol=0.0001, rtol=0.0001 ) - paddle.enable_static() if __name__ == '__main__': From 4ecda488eabf5ca4041d4d904f9f42cc04943c2f Mon Sep 17 00:00:00 2001 From: Micalling <2687920886@qq.com> Date: Wed, 9 Oct 2024 12:03:31 +0800 Subject: [PATCH 4/4] update --- test/legacy_test/test_get_window.py | 14 ++++++++++---- 1 file changed, 10 insertions(+), 4 deletions(-) diff --git a/test/legacy_test/test_get_window.py b/test/legacy_test/test_get_window.py index c793ade880821..189b45d257458 100644 --- a/test/legacy_test/test_get_window.py +++ b/test/legacy_test/test_get_window.py @@ -60,7 +60,6 @@ def test_window(self, window_type: str, n_fft: int): rtol=0.0001, ) - @parameterize([1, 512]) def test_window_and_exception(self, n_fft: int): window_scipy_gaussain = signal.windows.gaussian(n_fft, std=7) @@ -71,7 +70,10 @@ def test_window_and_exception(self, n_fft: int): window_scipy_gaussain, dtype=window_paddle_gaussian.dtype ) paddle.allclose( - window_scipy_gaussain, window_paddle_gaussian, atol=0.0001, rtol=0.0001 + window_scipy_gaussain, + window_paddle_gaussian, + atol=0.0001, + rtol=0.0001, ) window_scipy_general_gaussain = signal.windows.general_gaussian( @@ -81,10 +83,14 @@ def test_window_and_exception(self, n_fft: int): ('general_gaussian', 1, 7), n_fft, False ) window_scipy_general_gaussain = paddle.to_tensor( - window_scipy_general_gaussain, dtype=window_paddle_general_gaussian.dtype + window_scipy_general_gaussain, + dtype=window_paddle_general_gaussian.dtype, ) paddle.allclose( - window_scipy_gaussain, window_paddle_gaussian, atol=0.0001, rtol=0.0001 + window_scipy_gaussain, + window_paddle_gaussian, + atol=0.0001, + rtol=0.0001, ) window_scipy_exp = signal.windows.exponential(n_fft)