File size: 5,577 Bytes
434b0b0 | 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 | # Copyright (c) 2023-2024, Zexin He
#
# 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
#
# https://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 torch.nn as nn
from .modulate import ModLN
class BasicBlock(nn.Module):
"""
Transformer block that is in its simplest form.
Designed for PF-LRM architecture.
"""
# Block contains a self-attention layer and an MLP
def __init__(
self,
inner_dim: int,
num_heads: int,
eps: float,
attn_drop: float = 0.0,
attn_bias: bool = False,
mlp_ratio: float = 4.0,
mlp_drop: float = 0.0,
):
super().__init__()
self.norm1 = nn.LayerNorm(inner_dim, eps=eps)
self.self_attn = nn.MultiheadAttention(
embed_dim=inner_dim,
num_heads=num_heads,
dropout=attn_drop,
bias=attn_bias,
batch_first=True,
)
self.norm2 = nn.LayerNorm(inner_dim, eps=eps)
self.mlp = nn.Sequential(
nn.Linear(inner_dim, int(inner_dim * mlp_ratio)),
nn.GELU(),
nn.Dropout(mlp_drop),
nn.Linear(int(inner_dim * mlp_ratio), inner_dim),
nn.Dropout(mlp_drop),
)
def forward(self, x):
# x: [N, L, D]
before_sa = self.norm1(x)
x = x + self.self_attn(before_sa, before_sa, before_sa, need_weights=False)[0]
x = x + self.mlp(self.norm2(x))
return x
class ConditionBlock(nn.Module):
"""
Transformer block that takes in a cross-attention condition.
Designed for SparseLRM architecture.
"""
# Block contains a cross-attention layer, a self-attention layer, and an MLP
def __init__(
self,
inner_dim: int,
cond_dim: int,
num_heads: int,
eps: float,
attn_drop: float = 0.0,
attn_bias: bool = False,
mlp_ratio: float = 4.0,
mlp_drop: float = 0.0,
):
super().__init__()
self.norm1 = nn.LayerNorm(inner_dim, eps=eps)
self.cross_attn = nn.MultiheadAttention(
embed_dim=inner_dim,
num_heads=num_heads,
kdim=cond_dim,
vdim=cond_dim,
dropout=attn_drop,
bias=attn_bias,
batch_first=True,
)
self.norm2 = nn.LayerNorm(inner_dim, eps=eps)
self.self_attn = nn.MultiheadAttention(
embed_dim=inner_dim,
num_heads=num_heads,
dropout=attn_drop,
bias=attn_bias,
batch_first=True,
)
self.norm3 = nn.LayerNorm(inner_dim, eps=eps)
self.mlp = nn.Sequential(
nn.Linear(inner_dim, int(inner_dim * mlp_ratio)),
nn.GELU(),
nn.Dropout(mlp_drop),
nn.Linear(int(inner_dim * mlp_ratio), inner_dim),
nn.Dropout(mlp_drop),
)
def forward(self, x, cond):
# x: [N, L, D]
# cond: [N, L_cond, D_cond]
x = x + self.cross_attn(self.norm1(x), cond, cond, need_weights=False)[0]
before_sa = self.norm2(x)
x = x + self.self_attn(before_sa, before_sa, before_sa, need_weights=False)[0]
x = x + self.mlp(self.norm3(x))
return x
class ConditionModulationBlock(nn.Module):
"""
Transformer block that takes in a cross-attention condition and another modulation vector applied to sub-blocks.
Designed for raw LRM architecture.
"""
# Block contains a cross-attention layer, a self-attention layer, and an MLP
def __init__(
self,
inner_dim: int,
cond_dim: int,
mod_dim: int,
num_heads: int,
eps: float,
attn_drop: float = 0.0,
attn_bias: bool = False,
mlp_ratio: float = 4.0,
mlp_drop: float = 0.0,
):
super().__init__()
self.norm1 = ModLN(inner_dim, mod_dim, eps)
self.cross_attn = nn.MultiheadAttention(
embed_dim=inner_dim,
num_heads=num_heads,
kdim=cond_dim,
vdim=cond_dim,
dropout=attn_drop,
bias=attn_bias,
batch_first=True,
)
self.norm2 = ModLN(inner_dim, mod_dim, eps)
self.self_attn = nn.MultiheadAttention(
embed_dim=inner_dim,
num_heads=num_heads,
dropout=attn_drop,
bias=attn_bias,
batch_first=True,
)
self.norm3 = ModLN(inner_dim, mod_dim, eps)
self.mlp = nn.Sequential(
nn.Linear(inner_dim, int(inner_dim * mlp_ratio)),
nn.GELU(),
nn.Dropout(mlp_drop),
nn.Linear(int(inner_dim * mlp_ratio), inner_dim),
nn.Dropout(mlp_drop),
)
def forward(self, x, cond, mod):
# x: [N, L, D]
# cond: [N, L_cond, D_cond]
# mod: [N, D_mod]
x = x + self.cross_attn(self.norm1(x, mod), cond, cond, need_weights=False)[0]
before_sa = self.norm2(x, mod)
x = x + self.self_attn(before_sa, before_sa, before_sa, need_weights=False)[0]
x = x + self.mlp(self.norm3(x, mod))
return x
|