resnet_audio.py 13 KB
Newer Older
Sugon_ldc's avatar
Sugon_ldc committed
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
# Copyright (c) OpenMMLab. All rights reserved.
import torch.nn as nn
import torch.utils.checkpoint as cp
from mmcv.cnn import ConvModule, constant_init, kaiming_init
from mmcv.runner import load_checkpoint
from torch.nn.modules.batchnorm import _BatchNorm
from torch.nn.modules.utils import _ntuple

from ...utils import get_root_logger
from ..builder import BACKBONES


class Bottleneck2dAudio(nn.Module):
    """Bottleneck2D block for ResNet2D.

    Args:
        inplanes (int): Number of channels for the input in first conv3d layer.
        planes (int): Number of channels produced by some norm/conv3d layers.
        stride (int | tuple[int]): Stride in the conv layer. Default: 1.
        dilation (int): Spacing between kernel elements. Default: 1.
        downsample (nn.Module): Downsample layer. Default: None.
        factorize (bool): Whether to factorize kernel. Default: True.
        norm_cfg (dict):
            Config for norm layers. required keys are `type` and
            `requires_grad`. Default: None.
        with_cp (bool): Use checkpoint or not. Using checkpoint will save some
            memory while slowing down the training speed. Default: False.
    """
    expansion = 4

    def __init__(self,
                 inplanes,
                 planes,
                 stride=2,
                 dilation=1,
                 downsample=None,
                 factorize=True,
                 norm_cfg=None,
                 with_cp=False):
        super().__init__()

        self.inplanes = inplanes
        self.planes = planes
        self.stride = stride
        self.dilation = dilation
        self.factorize = factorize
        self.norm_cfg = norm_cfg
        self.with_cp = with_cp

        self.conv1_stride = 1
        self.conv2_stride = stride

        conv1_kernel_size = (1, 1)
        conv1_padding = 0
        conv2_kernel_size = (3, 3)
        conv2_padding = (dilation, dilation)
        self.conv1 = ConvModule(
            inplanes,
            planes,
            kernel_size=conv1_kernel_size,
            padding=conv1_padding,
            dilation=dilation,
            norm_cfg=self.norm_cfg,
            bias=False)
        self.conv2 = ConvModule(
            planes,
            planes,
            kernel_size=conv2_kernel_size,
            stride=stride,
            padding=conv2_padding,
            dilation=dilation,
            bias=False,
            conv_cfg=dict(type='ConvAudio') if factorize else dict(
                type='Conv'),
            norm_cfg=None,
            act_cfg=None)
        self.conv3 = ConvModule(
            2 * planes if factorize else planes,
            planes * self.expansion,
            kernel_size=1,
            bias=False,
            norm_cfg=self.norm_cfg,
            act_cfg=None)

        self.relu = nn.ReLU(inplace=True)
        self.downsample = downsample

    def forward(self, x):

        def _inner_forward(x):
            identity = x
            out = self.conv1(x)
            out = self.conv2(out)
            out = self.conv3(out)

            if self.downsample is not None:
                identity = self.downsample(x)
            out += identity

            return out

        if self.with_cp and x.requires_grad:
            out = cp.checkpoint(_inner_forward, x)
        else:
            out = _inner_forward(x)

        out = self.relu(out)

        return out


@BACKBONES.register_module()
class ResNetAudio(nn.Module):
    """ResNet 2d audio backbone. Reference:

        <https://arxiv.org/abs/2001.08740>`_.

    Args:
        depth (int): Depth of resnet, from {50, 101, 152}.
        pretrained (str | None): Name of pretrained model.
        in_channels (int): Channel num of input features. Default: 1.
        base_channels (int): Channel num of stem output features. Default: 32.
        num_stages (int): Resnet stages. Default: 4.
        strides (Sequence[int]): Strides of residual blocks of each stage.
            Default: (1, 2, 2, 2).
        dilations (Sequence[int]): Dilation of each stage.
            Default: (1, 1, 1, 1).
        conv1_kernel (int): Kernel size of the first conv layer. Default: 9.
        conv1_stride (int | tuple[int]): Stride of the first conv layer.
            Default: 1.
        frozen_stages (int): Stages to be frozen (all param fixed). -1 means
            not freezing any parameters.
        factorize (Sequence[int]): factorize Dims of each block for audio.
            Default: (1, 1, 0, 0).
        norm_eval (bool): Whether to set BN layers to eval mode, namely, freeze
            running stats (mean and var). Default: False.
        with_cp (bool): Use checkpoint or not. Using checkpoint will save some
            memory while slowing down the training speed. Default: False.
        conv_cfg (dict): Config for norm layers. Default: dict(type='Conv').
        norm_cfg (dict):
            Config for norm layers. required keys are `type` and
            `requires_grad`. Default: dict(type='BN2d', requires_grad=True).
        act_cfg (dict): Config for activate layers.
            Default: dict(type='ReLU', inplace=True).
        zero_init_residual (bool):
            Whether to use zero initialization for residual block,
            Default: True.
    """

    arch_settings = {
        # 18: (BasicBlock2dAudio, (2, 2, 2, 2)),
        # 34: (BasicBlock2dAudio, (3, 4, 6, 3)),
        50: (Bottleneck2dAudio, (3, 4, 6, 3)),
        101: (Bottleneck2dAudio, (3, 4, 23, 3)),
        152: (Bottleneck2dAudio, (3, 8, 36, 3))
    }

    def __init__(self,
                 depth,
                 pretrained,
                 in_channels=1,
                 num_stages=4,
                 base_channels=32,
                 strides=(1, 2, 2, 2),
                 dilations=(1, 1, 1, 1),
                 conv1_kernel=9,
                 conv1_stride=1,
                 frozen_stages=-1,
                 factorize=(1, 1, 0, 0),
                 norm_eval=False,
                 with_cp=False,
                 conv_cfg=dict(type='Conv'),
                 norm_cfg=dict(type='BN2d', requires_grad=True),
                 act_cfg=dict(type='ReLU', inplace=True),
                 zero_init_residual=True):
        super().__init__()
        if depth not in self.arch_settings:
            raise KeyError(f'invalid depth {depth} for resnet')
        self.depth = depth
        self.pretrained = pretrained
        self.in_channels = in_channels
        self.base_channels = base_channels
        self.num_stages = num_stages
        assert 1 <= num_stages <= 4
        self.dilations = dilations
        self.conv1_kernel = conv1_kernel
        self.conv1_stride = conv1_stride
        self.frozen_stages = frozen_stages
        self.stage_factorization = _ntuple(num_stages)(factorize)
        self.norm_eval = norm_eval
        self.with_cp = with_cp
        self.conv_cfg = conv_cfg
        self.norm_cfg = norm_cfg
        self.act_cfg = act_cfg
        self.zero_init_residual = zero_init_residual

        self.block, stage_blocks = self.arch_settings[depth]
        self.stage_blocks = stage_blocks[:num_stages]
        self.inplanes = self.base_channels

        self._make_stem_layer()

        self.res_layers = []
        for i, num_blocks in enumerate(self.stage_blocks):
            stride = strides[i]
            dilation = dilations[i]
            planes = self.base_channels * 2**i
            res_layer = self.make_res_layer(
                self.block,
                self.inplanes,
                planes,
                num_blocks,
                stride=stride,
                dilation=dilation,
                factorize=self.stage_factorization[i],
                norm_cfg=self.norm_cfg,
                with_cp=with_cp)
            self.inplanes = planes * self.block.expansion
            layer_name = f'layer{i + 1}'
            self.add_module(layer_name, res_layer)
            self.res_layers.append(layer_name)

        self.feat_dim = self.block.expansion * self.base_channels * 2**(
            len(self.stage_blocks) - 1)

    @staticmethod
    def make_res_layer(block,
                       inplanes,
                       planes,
                       blocks,
                       stride=1,
                       dilation=1,
                       factorize=1,
                       norm_cfg=None,
                       with_cp=False):
        """Build residual layer for ResNetAudio.

        Args:
            block (nn.Module): Residual module to be built.
            inplanes (int): Number of channels for the input feature
                in each block.
            planes (int): Number of channels for the output feature
                in each block.
            blocks (int): Number of residual blocks.
            stride (Sequence[int]): Strides of residual blocks of each stage.
                Default: (1, 2, 2, 2).
            dilation (int): Spacing between kernel elements. Default: 1.
            factorize (int | Sequence[int]): Determine whether to factorize
                for each block. Default: 1.
            norm_cfg (dict):
                Config for norm layers. required keys are `type` and
                `requires_grad`. Default: None.
            with_cp (bool): Use checkpoint or not. Using checkpoint will save
                some memory while slowing down the training speed.
                Default: False.

        Returns:
            A residual layer for the given config.
        """
        factorize = factorize if not isinstance(
            factorize, int) else (factorize, ) * blocks
        assert len(factorize) == blocks
        downsample = None
        if stride != 1 or inplanes != planes * block.expansion:
            downsample = ConvModule(
                inplanes,
                planes * block.expansion,
                kernel_size=1,
                stride=stride,
                bias=False,
                norm_cfg=norm_cfg,
                act_cfg=None)

        layers = []
        layers.append(
            block(
                inplanes,
                planes,
                stride,
                dilation,
                downsample,
                factorize=(factorize[0] == 1),
                norm_cfg=norm_cfg,
                with_cp=with_cp))
        inplanes = planes * block.expansion
        for i in range(1, blocks):
            layers.append(
                block(
                    inplanes,
                    planes,
                    1,
                    dilation,
                    factorize=(factorize[i] == 1),
                    norm_cfg=norm_cfg,
                    with_cp=with_cp))

        return nn.Sequential(*layers)

    def _make_stem_layer(self):
        """Construct the stem layers consists of a conv+norm+act module and a
        pooling layer."""
        self.conv1 = ConvModule(
            self.in_channels,
            self.base_channels,
            kernel_size=self.conv1_kernel,
            stride=self.conv1_stride,
            bias=False,
            conv_cfg=dict(type='ConvAudio', op='sum'),
            norm_cfg=self.norm_cfg,
            act_cfg=self.act_cfg)

    def _freeze_stages(self):
        """Prevent all the parameters from being optimized before
        ``self.frozen_stages``."""
        if self.frozen_stages >= 0:
            self.conv1.bn.eval()
            for m in [self.conv1.conv, self.conv1.bn]:
                for param in m.parameters():
                    param.requires_grad = False

        for i in range(1, self.frozen_stages + 1):
            m = getattr(self, f'layer{i}')
            m.eval()
            for param in m.parameters():
                param.requires_grad = False

    def init_weights(self):
        """Initiate the parameters either from existing checkpoint or from
        scratch."""
        if isinstance(self.pretrained, str):
            logger = get_root_logger()
            logger.info(f'load model from: {self.pretrained}')

            load_checkpoint(self, self.pretrained, strict=False, logger=logger)

        elif self.pretrained is None:
            for m in self.modules():
                if isinstance(m, nn.Conv2d):
                    kaiming_init(m)
                elif isinstance(m, _BatchNorm):
                    constant_init(m, 1)

            if self.zero_init_residual:
                for m in self.modules():
                    if isinstance(m, Bottleneck2dAudio):
                        constant_init(m.conv3.bn, 0)

        else:
            raise TypeError('pretrained must be a str or None')

    def forward(self, x):
        """Defines the computation performed at every call.

        Args:
            x (torch.Tensor): The input data.

        Returns:
            torch.Tensor: The feature of the input samples extracted
            by the backbone.
        """
        x = self.conv1(x)
        for layer_name in self.res_layers:
            res_layer = getattr(self, layer_name)
            x = res_layer(x)
        return x

    def train(self, mode=True):
        """Set the optimization status when training."""
        super().train(mode)
        self._freeze_stages()
        if mode and self.norm_eval:
            for m in self.modules():
                if isinstance(m, _BatchNorm):
                    m.eval()