Skip to content

Commit 3470da2

Browse files
rattus128adlerfaulkner
authored andcommitted
Reduce Peak WAN inference VRAM usage (comfyanonymous#9898)
* flux: Do the xq and xk ropes one at a time This was doing independendent interleaved tensor math on the q and k tensors, leading to the holding of more than the minimum intermediates in VRAM. On a bad day, it would VRAM OOM on xk intermediates. Do everything q and then everything k, so torch can garbage collect all of qs intermediates before k allocates its intermediates. This reduces peak VRAM usage for some WAN2.2 inferences (at least). * wan: Optimize qkv intermediates on attention As commented. The former logic computed independent pieces of QKV in parallel which help more inference intermediates in VRAM spiking VRAM usage. Fully roping Q and garbage collecting the intermediates before touching K reduces the peak inference VRAM usage.
1 parent 0814f13 commit 3470da2

File tree

2 files changed

+17
-14
lines changed

2 files changed

+17
-14
lines changed

comfy/ldm/flux/math.py

Lines changed: 5 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -35,11 +35,10 @@ def rope(pos: Tensor, dim: int, theta: int) -> Tensor:
3535
out = rearrange(out, "b n d (i j) -> b n d i j", i=2, j=2)
3636
return out.to(dtype=torch.float32, device=pos.device)
3737

38+
def apply_rope1(x: Tensor, freqs_cis: Tensor):
39+
x_ = x.to(dtype=freqs_cis.dtype).reshape(*x.shape[:-1], -1, 1, 2)
40+
x_out = freqs_cis[..., 0] * x_[..., 0] + freqs_cis[..., 1] * x_[..., 1]
41+
return x_out.reshape(*x.shape).type_as(x)
3842

3943
def apply_rope(xq: Tensor, xk: Tensor, freqs_cis: Tensor):
40-
xq_ = xq.to(dtype=freqs_cis.dtype).reshape(*xq.shape[:-1], -1, 1, 2)
41-
xk_ = xk.to(dtype=freqs_cis.dtype).reshape(*xk.shape[:-1], -1, 1, 2)
42-
xq_out = freqs_cis[..., 0] * xq_[..., 0] + freqs_cis[..., 1] * xq_[..., 1]
43-
xk_out = freqs_cis[..., 0] * xk_[..., 0] + freqs_cis[..., 1] * xk_[..., 1]
44-
return xq_out.reshape(*xq.shape).type_as(xq), xk_out.reshape(*xk.shape).type_as(xk)
45-
44+
return apply_rope1(xq, freqs_cis), apply_rope1(xk, freqs_cis)

comfy/ldm/wan/model.py

Lines changed: 12 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -8,7 +8,7 @@
88

99
from comfy.ldm.modules.attention import optimized_attention
1010
from comfy.ldm.flux.layers import EmbedND
11-
from comfy.ldm.flux.math import apply_rope
11+
from comfy.ldm.flux.math import apply_rope1
1212
import comfy.ldm.common_dit
1313
import comfy.model_management
1414
import comfy.patcher_extension
@@ -60,20 +60,24 @@ def forward(self, x, freqs, transformer_options={}):
6060
"""
6161
b, s, n, d = *x.shape[:2], self.num_heads, self.head_dim
6262

63-
# query, key, value function
64-
def qkv_fn(x):
63+
def qkv_fn_q(x):
6564
q = self.norm_q(self.q(x)).view(b, s, n, d)
65+
return apply_rope1(q, freqs)
66+
67+
def qkv_fn_k(x):
6668
k = self.norm_k(self.k(x)).view(b, s, n, d)
67-
v = self.v(x).view(b, s, n * d)
68-
return q, k, v
69+
return apply_rope1(k, freqs)
6970

70-
q, k, v = qkv_fn(x)
71-
q, k = apply_rope(q, k, freqs)
71+
#These two are VRAM hogs, so we want to do all of q computation and
72+
#have pytorch garbage collect the intermediates on the sub function
73+
#return before we touch k
74+
q = qkv_fn_q(x)
75+
k = qkv_fn_k(x)
7276

7377
x = optimized_attention(
7478
q.view(b, s, n * d),
7579
k.view(b, s, n * d),
76-
v,
80+
self.v(x).view(b, s, n * d),
7781
heads=self.num_heads,
7882
transformer_options=transformer_options,
7983
)

0 commit comments

Comments
 (0)