You can not select more than 25 topics Topics must start with a letter or number, can include dashes ('-') and can be up to 35 characters long.

modules.py 8.8KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257
  1. import torch
  2. from torch import nn
  3. from torch.nn import functional as F
  4. from einops import rearrange, repeat
  5. import math
  6. def scaled_dot_product(q, k, v, mask=None):
  7. d_k = q.size()[-1]
  8. attn_logits = torch.matmul(q, k.transpose(-2, -1))
  9. attn_logits = attn_logits / math.sqrt(d_k)
  10. if mask is not None:
  11. attn_logits = attn_logits.masked_fill(mask == 0, -9e15)
  12. attention = F.softmax(attn_logits, dim=-1)
  13. values = torch.matmul(attention, v)
  14. return values, attention
  15. """ Fusion Module"""
  16. class ASM(nn.Module):
  17. def __init__(self, in_channels, all_channels):
  18. super(ASM, self).__init__()
  19. self.non_local = NonLocalBlock(in_channels)
  20. def forward(self, lc, fuse, gc):
  21. fuse = self.non_local(fuse)
  22. fuse = torch.cat([lc, fuse, gc], dim=1)
  23. return fuse
  24. """
  25. Squeeze and Excitation Layer
  26. https://arxiv.org/abs/1709.01507
  27. """
  28. class SELayer(nn.Module):
  29. def __init__(self, channel, reduction=16):
  30. super(SELayer, self).__init__()
  31. self.avg_pool = nn.AdaptiveAvgPool2d(1)
  32. self.fc = nn.Sequential(
  33. nn.Linear(channel, channel // reduction, bias=False),
  34. nn.ReLU(inplace=True),
  35. nn.Linear(channel // reduction, channel, bias=False),
  36. nn.Sigmoid()
  37. )
  38. def forward(self, x):
  39. b, c, _, _ = x.size()
  40. y = self.avg_pool(x).view(b, c)
  41. y = self.fc(y).view(b, c, 1, 1)
  42. return x * y.expand_as(x)
  43. """
  44. Non Local Block
  45. https://arxiv.org/abs/1711.07971
  46. """
  47. class NonLocalBlock(nn.Module):
  48. def __init__(self, in_channels, inter_channels=None, sub_sample=True, bn_layer=True):
  49. super(NonLocalBlock, self).__init__()
  50. self.sub_sample = sub_sample
  51. self.in_channels = in_channels
  52. self.inter_channels = inter_channels
  53. if self.inter_channels is None:
  54. self.inter_channels = in_channels // 2
  55. if self.inter_channels == 0:
  56. self.inter_channels = 1
  57. self.g = nn.Conv2d(in_channels=self.in_channels, out_channels=self.inter_channels,
  58. kernel_size=1, stride=1, padding=0)
  59. if bn_layer:
  60. self.W = nn.Sequential(
  61. nn.Conv2d(in_channels=self.inter_channels, out_channels=self.in_channels,
  62. kernel_size=1, stride=1, padding=0),
  63. nn.BatchNorm2d(self.in_channels)
  64. )
  65. nn.init.constant_(self.W[1].weight, 0)
  66. nn.init.constant_(self.W[1].bias, 0)
  67. else:
  68. self.W = nn.Conv2d(in_channels=self.inter_channels, out_channels=self.in_channels,
  69. kernel_size=1, stride=1, padding=0)
  70. nn.init.constant_(self.W.weight, 0)
  71. nn.init.constant_(self.W.bias, 0)
  72. self.theta = nn.Conv2d(in_channels=self.in_channels, out_channels=self.inter_channels,
  73. kernel_size=1, stride=1, padding=0)
  74. self.phi = nn.Conv2d(in_channels=self.in_channels, out_channels=self.inter_channels,
  75. kernel_size=1, stride=1, padding=0)
  76. if sub_sample:
  77. self.g = nn.Sequential(self.g, nn.MaxPool2d(kernel_size=(2, 2)))
  78. self.phi = nn.Sequential(self.phi, nn.MaxPool2d(kernel_size=(2, 2)))
  79. def forward(self, x):
  80. batch_size = x.size(0)
  81. g_x = self.g(x).view(batch_size, self.inter_channels, -1)
  82. g_x = g_x.permute(0, 2, 1)
  83. theta_x = self.theta(x).view(batch_size, self.inter_channels, -1)
  84. theta_x = theta_x.permute(0, 2, 1)
  85. phi_x = self.phi(x).view(batch_size, self.inter_channels, -1)
  86. f = torch.matmul(theta_x, phi_x)
  87. f_div_C = F.softmax(f, dim=-1)
  88. y = torch.matmul(f_div_C, g_x)
  89. y = y.permute(0, 2, 1).contiguous()
  90. y = y.view(batch_size, self.inter_channels, *x.size()[2:])
  91. W_y = self.W(y)
  92. z = W_y + x
  93. return z
  94. #AGCM Module
  95. class CrossNonLocalBlock(nn.Module):
  96. def __init__(self, in_channels_source,in_channels_target, inter_channels, sub_sample=False, bn_layer=True):
  97. super(CrossNonLocalBlock, self).__init__()
  98. self.sub_sample = sub_sample
  99. self.in_channels_source = in_channels_source
  100. self.in_channels_target = in_channels_target
  101. self.inter_channels = inter_channels
  102. """
  103. if self.inter_channels is None:
  104. self.inter_channels = in_channels // 2
  105. if self.inter_channels == 0:
  106. self.inter_channels = 1
  107. """
  108. self.g = nn.Conv2d(in_channels=self.in_channels_source, out_channels=self.inter_channels,
  109. kernel_size=1, stride=1, padding=0)
  110. self.theta = nn.Conv2d(in_channels=self.in_channels_source, out_channels=self.inter_channels,
  111. kernel_size=1, stride=1, padding=0)
  112. self.phi = nn.Conv2d(in_channels=self.in_channels_target, out_channels=self.inter_channels,
  113. kernel_size=1, stride=1, padding=0)
  114. if bn_layer:
  115. self.W = nn.Sequential(
  116. nn.Conv2d(in_channels=self.inter_channels, out_channels=self.in_channels_target,
  117. kernel_size=1, stride=1, padding=0),
  118. nn.BatchNorm2d(self.in_channels_target)
  119. )
  120. nn.init.constant_(self.W[1].weight, 0)
  121. nn.init.constant_(self.W[1].bias, 0)
  122. if sub_sample:
  123. self.g = nn.Sequential(self.g, nn.MaxPool2d(kernel_size=(2, 2)))
  124. self.phi = nn.Sequential(self.phi, nn.MaxPool2d(kernel_size=(2, 2)))
  125. def forward(self,x,l):
  126. batch_size = x.size(0)
  127. g_x = self.g(x).view(batch_size, self.inter_channels, -1)
  128. g_x = g_x.permute(0, 2, 1) #source
  129. theta_x1 = self.theta(x)
  130. theta_x = self.theta(x).view(batch_size, self.inter_channels, -1)
  131. theta_x = theta_x.permute(0, 2, 1) #source
  132. phi_x = self.phi(l).view(batch_size, self.inter_channels, -1) #target
  133. f = torch.matmul(theta_x, phi_x)
  134. f_div_C = F.softmax(f, dim=-1)
  135. f_div_C = f_div_C.permute(0,2,1)
  136. y = torch.matmul(f_div_C, g_x)
  137. y = y.permute(0, 2, 1).contiguous()
  138. y = y.view(batch_size, self.inter_channels, *l.size()[2:])
  139. W_y = self.W(y)
  140. z = W_y + l
  141. return z
  142. #SFEM module
  143. class NonLocalBlock_PatchWise(nn.Module):
  144. def __init__(self, in_channel, inter_channel, patch_factor):
  145. super(NonLocalBlock_PatchWise, self).__init__()
  146. "Embedding dimension must be 0 modulo number of heads."
  147. self.in_channel = in_channel
  148. self.patch_factor = patch_factor
  149. self.patch_width = int(8/self.patch_factor)
  150. self.patch_height = int(8/self.patch_factor)
  151. self.stride_width = int(8/self.patch_factor)
  152. self.stride_height = int(8/self.patch_factor)
  153. self.unfold = nn.Unfold(kernel_size=(self.patch_width, self.patch_height), stride=(self.stride_width, self.stride_height))
  154. self.adp = nn.AdaptiveAvgPool2d(8)
  155. self.bottleneck = nn.Conv2d(64,inter_channel,kernel_size=(1,1))
  156. self.non_block = NonLocalBlock(self.in_channel)
  157. self.adp_post = nn.AdaptiveAvgPool2d((8,8))
  158. def forward(self, x):
  159. batch_size = x.size(0)
  160. x_up = self.adp(x)
  161. x_up = self.unfold(x)
  162. batch_size,p_dim,p_size = x_up.size()
  163. x_up = x_up.view(batch_size,-1,self.in_channel,p_size)
  164. final_output = torch.tensor([]).cuda()
  165. index = torch.range(0,p_size,1,dtype=torch.int64).cuda()
  166. for i in range(int(p_size)):
  167. divide = torch.index_select(x_up, 3, index[i])
  168. divide = divide.view(batch_size,-1,self.in_channel)
  169. patch_width = int(divide.size(1) ** 0.5)
  170. divide = divide.reshape(batch_size,self.in_channel,patch_width,patch_width) # tensor to operate on
  171. attn = self.non_block(divide)
  172. output = attn.view(batch_size,-1,self.in_channel,1)
  173. final_output = torch.cat((final_output,output),dim=3)
  174. final_output = final_output.view(batch_size, self.in_channel, 8,8)
  175. return final_output
  176. class GCM_up(nn.Module):
  177. def __init__(self, in_channels, out_channels):
  178. super(GCM_up, self).__init__()
  179. self.adp = nn.AdaptiveAvgPool2d((8,8))
  180. self.patch1 = NonLocalBlock_PatchWise(in_channels,out_channels,2)
  181. self.patch2 = NonLocalBlock_PatchWise(in_channels,out_channels,4)
  182. self.patch3 = NonLocalBlock(256,64)
  183. self.fuse = SELayer(3*256)
  184. self.conv = nn.Conv2d(3*256, out_channels, 1, 1)
  185. self.relu = nn.ReLU(inplace=True)
  186. def forward(self, x):
  187. b,c,h,w = x.size()
  188. x = self.adp(x)
  189. patch1 = self.patch1(x)
  190. patch2 = self.patch2(x)
  191. patch3 = self.patch3(x)
  192. global_cat = torch.cat((patch1, patch2, patch3), dim=1)
  193. fuse = self.relu(self.conv(self.fuse(global_cat)))
  194. adp_post = nn.AdaptiveAvgPool2d((h,w))
  195. fuse = adp_post(fuse)
  196. return fuse