dilated_encoder.py 3.9 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109
  1. # Copyright (c) OpenMMLab. All rights reserved.
  2. import torch.nn as nn
  3. from mmcv.cnn import ConvModule, is_norm
  4. from mmengine.model import caffe2_xavier_init, constant_init, normal_init
  5. from torch.nn import BatchNorm2d
  6. from mmdet.registry import MODELS
  7. class Bottleneck(nn.Module):
  8. """Bottleneck block for DilatedEncoder used in `YOLOF.
  9. <https://arxiv.org/abs/2103.09460>`.
  10. The Bottleneck contains three ConvLayers and one residual connection.
  11. Args:
  12. in_channels (int): The number of input channels.
  13. mid_channels (int): The number of middle output channels.
  14. dilation (int): Dilation rate.
  15. norm_cfg (dict): Dictionary to construct and config norm layer.
  16. """
  17. def __init__(self,
  18. in_channels,
  19. mid_channels,
  20. dilation,
  21. norm_cfg=dict(type='BN', requires_grad=True)):
  22. super(Bottleneck, self).__init__()
  23. self.conv1 = ConvModule(
  24. in_channels, mid_channels, 1, norm_cfg=norm_cfg)
  25. self.conv2 = ConvModule(
  26. mid_channels,
  27. mid_channels,
  28. 3,
  29. padding=dilation,
  30. dilation=dilation,
  31. norm_cfg=norm_cfg)
  32. self.conv3 = ConvModule(
  33. mid_channels, in_channels, 1, norm_cfg=norm_cfg)
  34. def forward(self, x):
  35. identity = x
  36. out = self.conv1(x)
  37. out = self.conv2(out)
  38. out = self.conv3(out)
  39. out = out + identity
  40. return out
  41. @MODELS.register_module()
  42. class DilatedEncoder(nn.Module):
  43. """Dilated Encoder for YOLOF <https://arxiv.org/abs/2103.09460>`.
  44. This module contains two types of components:
  45. - the original FPN lateral convolution layer and fpn convolution layer,
  46. which are 1x1 conv + 3x3 conv
  47. - the dilated residual block
  48. Args:
  49. in_channels (int): The number of input channels.
  50. out_channels (int): The number of output channels.
  51. block_mid_channels (int): The number of middle block output channels
  52. num_residual_blocks (int): The number of residual blocks.
  53. block_dilations (list): The list of residual blocks dilation.
  54. """
  55. def __init__(self, in_channels, out_channels, block_mid_channels,
  56. num_residual_blocks, block_dilations):
  57. super(DilatedEncoder, self).__init__()
  58. self.in_channels = in_channels
  59. self.out_channels = out_channels
  60. self.block_mid_channels = block_mid_channels
  61. self.num_residual_blocks = num_residual_blocks
  62. self.block_dilations = block_dilations
  63. self._init_layers()
  64. def _init_layers(self):
  65. self.lateral_conv = nn.Conv2d(
  66. self.in_channels, self.out_channels, kernel_size=1)
  67. self.lateral_norm = BatchNorm2d(self.out_channels)
  68. self.fpn_conv = nn.Conv2d(
  69. self.out_channels, self.out_channels, kernel_size=3, padding=1)
  70. self.fpn_norm = BatchNorm2d(self.out_channels)
  71. encoder_blocks = []
  72. for i in range(self.num_residual_blocks):
  73. dilation = self.block_dilations[i]
  74. encoder_blocks.append(
  75. Bottleneck(
  76. self.out_channels,
  77. self.block_mid_channels,
  78. dilation=dilation))
  79. self.dilated_encoder_blocks = nn.Sequential(*encoder_blocks)
  80. def init_weights(self):
  81. caffe2_xavier_init(self.lateral_conv)
  82. caffe2_xavier_init(self.fpn_conv)
  83. for m in [self.lateral_norm, self.fpn_norm]:
  84. constant_init(m, 1)
  85. for m in self.dilated_encoder_blocks.modules():
  86. if isinstance(m, nn.Conv2d):
  87. normal_init(m, mean=0, std=0.01)
  88. if is_norm(m):
  89. constant_init(m, 1)
  90. def forward(self, feature):
  91. out = self.lateral_norm(self.lateral_conv(feature[-1]))
  92. out = self.fpn_norm(self.fpn_conv(out))
  93. return self.dilated_encoder_blocks(out),