Skip to content
Merged
Show file tree
Hide file tree
Changes from all commits
Commits
Show all changes
61 commits
Select commit Hold shift + click to select a range
dceb05c
initial commit
NouamaneTazi Aug 29, 2022
e7ff942
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Aug 29, 2022
45696fd
make UNet stream capturable
NouamaneTazi Aug 29, 2022
d5edb19
try to fix noise_pred value
NouamaneTazi Aug 29, 2022
89143b1
remove cuda graph and keep NB
NouamaneTazi Aug 30, 2022
2f4a34b
non blocking unet with PNDMScheduler
NouamaneTazi Aug 30, 2022
e6e41ae
make timesteps np arrays for pndm scheduler
NouamaneTazi Aug 30, 2022
95051ae
make max async in pndm
NouamaneTazi Aug 30, 2022
0b85b6f
use channel last format in unet
NouamaneTazi Sep 6, 2022
101b8b0
avoid moving timesteps device in each unet call
NouamaneTazi Sep 6, 2022
493a64a
avoid memcpy op in `get_timestep_embedding`
NouamaneTazi Sep 6, 2022
75b4f0c
add `channels_last` kwarg to `DiffusionPipeline.from_pretrained`
NouamaneTazi Sep 6, 2022
1639f69
update TODO
NouamaneTazi Sep 6, 2022
f2176e9
replace `channels_last` kwarg with `memory_format` for more generality
NouamaneTazi Sep 6, 2022
08db0c3
revert the channels_last changes to leave it for another PR
NouamaneTazi Sep 8, 2022
de714e7
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Sep 8, 2022
98a5301
Merge branch 'main' of https://github.com/NouamaneTazi/diffusers into…
NouamaneTazi Sep 13, 2022
9647752
remove non_blocking when moving input ids to device
NouamaneTazi Sep 13, 2022
cab7b28
remove blocking from all .to() operations at beginning of pipeline
NouamaneTazi Sep 13, 2022
acb8397
fix merging
NouamaneTazi Sep 13, 2022
39994cc
fix merging
NouamaneTazi Sep 13, 2022
d30f968
model can run in other precisions without autocast
NouamaneTazi Sep 13, 2022
0c70c0e
attn refactoring
NouamaneTazi Sep 14, 2022
e422eb3
Revert "attn refactoring"
NouamaneTazi Sep 14, 2022
4e67675
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Sep 19, 2022
cec5928
remove restriction to run conv_norm in fp32
NouamaneTazi Sep 21, 2022
c0dd0e9
use `baddbmm` instead of `matmul`for better in attention for better perf
NouamaneTazi Sep 21, 2022
006ccb8
removing all reshapes to test perf
NouamaneTazi Sep 21, 2022
75fa029
Revert "removing all reshapes to test perf"
NouamaneTazi Sep 21, 2022
31c58ea
add shapes comments
NouamaneTazi Sep 21, 2022
2fa9c69
hardcore whats needed for jitting
NouamaneTazi Sep 22, 2022
47c668c
Revert "hardcore whats needed for jitting"
NouamaneTazi Sep 22, 2022
cc9bc13
Revert "remove restriction to run conv_norm in fp32"
NouamaneTazi Sep 22, 2022
419fde3
revert using baddmm in attention's forward
NouamaneTazi Sep 22, 2022
9312809
cleanup comment
NouamaneTazi Sep 22, 2022
03a2ee7
remove restriction to run conv_norm in fp32. no quality loss was noticed
NouamaneTazi Sep 22, 2022
d0b5579
add more optimizations techniques to docs
NouamaneTazi Sep 22, 2022
3bdf1ed
Revert "add shapes comments"
NouamaneTazi Sep 22, 2022
aeddb45
apply suggestions
NouamaneTazi Sep 24, 2022
7d0347f
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Sep 24, 2022
f409172
make quality
NouamaneTazi Sep 24, 2022
76dda3e
apply suggestions
NouamaneTazi Sep 27, 2022
8929d76
styling
NouamaneTazi Sep 27, 2022
318bdd0
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Sep 27, 2022
98e80da
`scheduler.timesteps` are now arrays so we dont need .to()
NouamaneTazi Sep 27, 2022
9b1ec08
remove useless .type()
NouamaneTazi Sep 28, 2022
61ed4c3
use mean instead of max in `test_stable_diffusion_inpaint_pipeline_k_…
NouamaneTazi Sep 28, 2022
fe6eed5
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Sep 28, 2022
e3c38e8
move scheduler timestamps to correct device if tensors
NouamaneTazi Sep 28, 2022
f25f1c1
add device to `set_timesteps` in LMSD scheduler
NouamaneTazi Sep 29, 2022
00d5a51
`self.scheduler.set_timesteps` now uses device arg for schedulers tha…
NouamaneTazi Sep 29, 2022
0cd4613
quick fix
NouamaneTazi Sep 29, 2022
0fb42d4
styling
NouamaneTazi Sep 29, 2022
e6969ed
remove kwargs from schedulers `set_timesteps`
NouamaneTazi Sep 29, 2022
e011d15
Merge branch 'main' of https://github.com/huggingface/diffusers into …
NouamaneTazi Sep 29, 2022
2ad3353
revert to using max in K-LMS inpaint pipeline test
NouamaneTazi Sep 29, 2022
7183202
Revert "`self.scheduler.set_timesteps` now uses device arg for schedu…
NouamaneTazi Sep 29, 2022
da67fe6
move timesteps to correct device before loop in SD pipeline
NouamaneTazi Sep 29, 2022
c8cc2ba
apply previous fix to other SD pipelines
NouamaneTazi Sep 29, 2022
b6162dc
UNet now accepts tensor timesteps even on wrong device, to avoid errors
NouamaneTazi Sep 29, 2022
9a1fb03
fix pipeline when timesteps are arrays with strides
NouamaneTazi Sep 29, 2022
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
198 changes: 195 additions & 3 deletions docs/source/optimization/fp16.mdx
Original file line number Diff line number Diff line change
Expand Up @@ -14,7 +14,64 @@ specific language governing permissions and limitations under the License.

We present some techniques and ideas to optimize 🤗 Diffusers _inference_ for memory or speed.

## CUDA `autocast`
<table>
<tr>
<td>
<td>Latency
<td>Speedup
<tr>
<tr>
<td>original
<td>9.50s
<td>x1
<tr>
<tr>
<td>cuDNN auto-tuner
<td>9.37s
<td>x1.01
<tr>
<td>autocast (fp16)
<td>5.47s
<td>x1.91
<tr>
<td>fp16
<td>3.61s
<td>x2.91
<tr>
<td>channels last
<td>3.30s
<td>x2.87
<tr>
<tr>
<td>traced UNet
<td>3.21s
<td>x2.96
</table>
<em>obtained on NVIDIA TITAN RTX by generating a single image of size 512x512 from the prompt "a photo of an astronaut riding a horse on mars" with 50 DDIM steps.</em>

## Enable cuDNN auto-tuner

[NVIDIA cuDNN](https://developer.nvidia.com/cudnn) supports many algorithms to compute a convolution. Autotuner runs a short benchmark and selects the kernel with the best performance on a given hardware for a given input size.

Since we’re using **convolutional networks** (other types currently not supported), we can enable cuDNN autotuner before launching the inference by setting:

```python
import torch

torch.backends.cudnn.benchmark = True
```

### Use tf32 instead of fp32 (on Ampere and later CUDA devices)

On Ampere and later CUDA devices matrix multiplications and convolutions can use the TensorFloat32 (TF32) mode for faster but slightly less accurate computations. By default PyTorch enables TF32 mode for convolutions but not matrix multiplications, and unless a network requires full float32 precision we recommend enabling this setting for matrix multiplications, too. It can significantly speed up computations with typically negligible loss of numerical accuracy. You can read more about it [here](https://huggingface.co/docs/transformers/v4.18.0/en/performance#tf32). All you need to do is to add this before your inference:

```python
import torch

torch.backends.cuda.matmul.allow_tf32 = True
```

## Automatic mixed precision (AMP)

If you use a CUDA GPU, you can take advantage of `torch.autocast` to perform inference roughly twice as fast at the cost of slightly lower precision. All you need to do is put your inference call inside an `autocast` context manager. The following example shows how to do it using Stable Diffusion text-to-image generation as an example:

Expand Down Expand Up @@ -47,7 +104,7 @@ pipe = StableDiffusionPipeline.from_pretrained(

## Sliced attention for additional memory savings

For even additional memory savings, you can use a sliced version of attention that performs the computation in steps instead of all at once.
For even additional memory savings, you can use a sliced version of attention that performs the computation in steps instead of all at once.

<Tip>
Attention slicing is useful even if a batch size of just 1 is used - as long as the model uses more than one attention head. If there is more than one attention head the *QK^T* attention matrix can be computed sequentially for each head which can save a significant amount of memory.
Expand All @@ -73,4 +130,139 @@ with torch.autocast("cuda"):
image = pipe(prompt).images[0]
```

There's a small performance penalty of about 10% slower inference times, but this method allows you to use Stable Diffusion in as little as 3.2 GB of VRAM!
There's a small performance penalty of about 10% slower inference times, but this method allows you to use Stable Diffusion in as little as 3.2 GB of VRAM!

## Using Channels Last memory format

Channels last memory format is an alternative way of ordering NCHW tensors in memory preserving dimensions ordering. Channels last tensors ordered in such a way that channels become the densest dimension (aka storing images pixel-per-pixel). Since not all operators currently support channels last format it may result in a worst performance, so it's better to try it and see if it works for your model.

For example, in order to set the UNet model in our pipeline to use channels last format, we can use the following:

```python
print(pipe.unet.conv_out.state_dict()["weight"].stride()) # (2880, 9, 3, 1)
pipe.unet.to(memory_format=torch.channels_last) # in-place operation
print(
pipe.unet.conv_out.state_dict()["weight"].stride()
) # (2880, 1, 960, 320) haveing a stride of 1 for the 2nd dimension proves that it works
```

## Tracing

Tracing runs an example input tensor through your model, and captures the operations that are invoked as that input makes its way through the model's layers so that an executable or `ScriptFunction` is returned that will be optimized using just-in-time compilation.

To trace our UNet model, we can use the following:

```python
import time
import torch
from diffusers import StableDiffusionPipeline
import functools

# torch disable grad
torch.set_grad_enabled(False)

# set variables
n_experiments = 2
unet_runs_per_experiment = 50

# load inputs
def generate_inputs():
sample = torch.randn(2, 4, 64, 64).half().cuda()
timestep = torch.rand(1).half().cuda() * 999
encoder_hidden_states = torch.randn(2, 77, 768).half().cuda()
return sample, timestep, encoder_hidden_states


pipe = StableDiffusionPipeline.from_pretrained(
"CompVis/stable-diffusion-v1-4",
# scheduler=scheduler,
use_auth_token=True,
revision="fp16",
torch_dtype=torch.float16,
).to("cuda")
unet = pipe.unet
unet.eval()
unet.to(memory_format=torch.channels_last) # use channels_last memory format
unet.forward = functools.partial(unet.forward, return_dict=False) # set return_dict=False as default

# warmup
for _ in range(3):
with torch.inference_mode():
inputs = generate_inputs()
orig_output = unet(*inputs)

# trace
print("tracing..")
unet_traced = torch.jit.trace(unet, inputs)
unet_traced.eval()
print("done tracing")


# warmup and optimize graph
for _ in range(5):
with torch.inference_mode():
inputs = generate_inputs()
orig_output = unet_traced(*inputs)


# benchmarking
with torch.inference_mode():
for _ in range(n_experiments):
torch.cuda.synchronize()
start_time = time.time()
for _ in range(unet_runs_per_experiment):
orig_output = unet_traced(*inputs)
torch.cuda.synchronize()
print(f"unet traced inference took {time.time() - start_time:.2f} seconds")
for _ in range(n_experiments):
torch.cuda.synchronize()
start_time = time.time()
for _ in range(unet_runs_per_experiment):
orig_output = unet(*inputs)
torch.cuda.synchronize()
print(f"unet inference took {time.time() - start_time:.2f} seconds")

# save the model
unet_traced.save("unet_traced.pt")
```

Then we can replace the `unet` attribute of the pipeline with the traced model like the following

```python
from diffusers import StableDiffusionPipeline
import torch
from dataclasses import dataclass


@dataclass
class UNet2DConditionOutput:
sample: torch.FloatTensor


pipe = StableDiffusionPipeline.from_pretrained(
"CompVis/stable-diffusion-v1-4",
# scheduler=scheduler,
use_auth_token=True,
revision="fp16",
torch_dtype=torch.float16,
).to("cuda")

# use jitted unet
unet_traced = torch.jit.load("unet_traced.pt")
# del pipe.unet
class TracedUNet(torch.nn.Module):
def __init__(self):
super().__init__()
self.in_channels = pipe.unet.in_channels
self.device = pipe.unet.device

def forward(self, latent_model_input, t, encoder_hidden_states):
sample = unet_traced(latent_model_input, t, encoder_hidden_states)[0]
return UNet2DConditionOutput(sample=sample)


pipe.unet = TracedUNet()

with torch.inference_mode():
image = pipe([prompt] * 1, num_inference_steps=50).images[0]
```
15 changes: 11 additions & 4 deletions src/diffusers/models/attention.py
Original file line number Diff line number Diff line change
Expand Up @@ -72,8 +72,7 @@ def forward(self, hidden_states):

# get scores
scale = 1 / math.sqrt(math.sqrt(self.channels / self.num_heads))

attention_scores = torch.matmul(query_states * scale, key_states.transpose(-1, -2) * scale)
attention_scores = torch.matmul(query_states * scale, key_states.transpose(-1, -2) * scale) # TODO: use baddmm
attention_probs = torch.softmax(attention_scores.float(), dim=-1).type(attention_scores.dtype)

# compute attention output
Expand Down Expand Up @@ -275,7 +274,13 @@ def forward(self, hidden_states, context=None, mask=None):
return self.to_out(hidden_states)

def _attention(self, query, key, value):
attention_scores = torch.matmul(query, key.transpose(-1, -2)) * self.scale
attention_scores = torch.baddbmm(
torch.empty(query.shape[0], query.shape[1], key.shape[1], dtype=query.dtype, device=query.device),
query,
key.transpose(-1, -2),
beta=0,
alpha=self.scale,
)
attention_probs = attention_scores.softmax(dim=-1)
# compute attention output
hidden_states = torch.matmul(attention_probs, value)
Expand All @@ -292,7 +297,9 @@ def _sliced_attention(self, query, key, value, sequence_length, dim):
for i in range(hidden_states.shape[0] // slice_size):
start_idx = i * slice_size
end_idx = (i + 1) * slice_size
attn_slice = torch.matmul(query[start_idx:end_idx], key[start_idx:end_idx].transpose(1, 2)) * self.scale
attn_slice = (
torch.matmul(query[start_idx:end_idx], key[start_idx:end_idx].transpose(1, 2)) * self.scale
) # TODO: use baddbmm for better performance
attn_slice = attn_slice.softmax(dim=-1)
attn_slice = torch.matmul(attn_slice, value[start_idx:end_idx])

Expand Down
6 changes: 4 additions & 2 deletions src/diffusers/models/embeddings.py
Original file line number Diff line number Diff line change
Expand Up @@ -37,10 +37,12 @@ def get_timestep_embedding(
assert len(timesteps.shape) == 1, "Timesteps should be a 1d-array"

half_dim = embedding_dim // 2
exponent = -math.log(max_period) * torch.arange(start=0, end=half_dim, dtype=torch.float32)
exponent = -math.log(max_period) * torch.arange(
start=0, end=half_dim, dtype=torch.float32, device=timesteps.device
Comment on lines +40 to +41
Copy link
Member Author

@NouamaneTazi NouamaneTazi Sep 29, 2022

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Do we need this opereration to run in fp32 even when the pipeline runs in fp16?

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Happy to try it out - let's maybe do it in a follow-up PR? :-)

)
exponent = exponent / (half_dim - downscale_freq_shift)

emb = torch.exp(exponent).to(device=timesteps.device)
emb = torch.exp(exponent)
emb = timesteps[:, None].float() * emb[None, :]
Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Same as previous comment


# scale embeddings
Expand Down
4 changes: 2 additions & 2 deletions src/diffusers/models/resnet.py
Original file line number Diff line number Diff line change
Expand Up @@ -331,7 +331,7 @@ def forward(self, x, temb):

# make sure hidden states is in float32
# when running in half-precision
hidden_states = self.norm1(hidden_states).type(hidden_states.dtype)
hidden_states = self.norm1(hidden_states)
hidden_states = self.nonlinearity(hidden_states)

if self.upsample is not None:
Expand All @@ -349,7 +349,7 @@ def forward(self, x, temb):

# make sure hidden states is in float32
# when running in half-precision
hidden_states = self.norm2(hidden_states).type(hidden_states.dtype)
hidden_states = self.norm2(hidden_states)
hidden_states = self.nonlinearity(hidden_states)

hidden_states = self.dropout(hidden_states)
Expand Down
8 changes: 4 additions & 4 deletions src/diffusers/models/unet_2d_condition.py
Original file line number Diff line number Diff line change
Expand Up @@ -230,16 +230,16 @@ def forward(
# 1. time
timesteps = timestep
if not torch.is_tensor(timesteps):
# TODO: this requires sync between CPU and GPU. So try to pass timesteps as tensors if you can
timesteps = torch.tensor([timesteps], dtype=torch.long, device=sample.device)
Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Do we need int64 tensors for timesteps, no matter the pipeline's precision?

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

At least for all the Stable Diffusion applications I've seen so far, timesteps are ints in the range 0..1000.

Even if other diffusion models do several orders of magnitude more than that, you'd think torch.int would be plenty.

Unless there's some byte alignment optimization reason to specifically make them 64-bit?

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

I think we can leave timesteps as int32

elif torch.is_tensor(timesteps) and len(timesteps.shape) == 0:
timesteps = timesteps.to(dtype=torch.float32)
timesteps = timesteps[None].to(device=sample.device)
timesteps = timesteps[None].to(sample.device)

# broadcast to batch dimension in a way that's compatible with ONNX/Core ML
timesteps = timesteps.expand(sample.shape[0])

t_emb = self.time_proj(timesteps)
emb = self.time_embedding(t_emb)
emb = self.time_embedding(t_emb.to(self.dtype))

# 2. pre-process
sample = self.conv_in(sample)
Expand Down Expand Up @@ -279,7 +279,7 @@ def forward(
# 6. post-process
# make sure hidden states is in float32
# when running in half-precision
sample = self.conv_norm_out(sample.float()).type(sample.dtype)
sample = self.conv_norm_out(sample)
sample = self.conv_act(sample)
sample = self.conv_out(sample)

Expand Down
Original file line number Diff line number Diff line change
Expand Up @@ -225,15 +225,23 @@ def __call__(
latents_shape,
generator=generator,
device=latents_device,
dtype=text_embeddings.dtype,
)
else:
if latents.shape != latents_shape:
raise ValueError(f"Unexpected latents shape, got {latents.shape}, expected {latents_shape}")
latents = latents.to(self.device)
latents = latents.to(latents_device)

# set timesteps
self.scheduler.set_timesteps(num_inference_steps)

# Some schedulers like PNDM have timesteps as arrays
# It's more optimzed to move all timesteps to correct device beforehand
if torch.is_tensor(self.scheduler.timesteps):
timesteps_tensor = self.scheduler.timesteps.to(self.device)
else:
timesteps_tensor = torch.tensor(self.scheduler.timesteps.copy(), device=self.device)
Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Should we make the dtype for timesteps int64?


# if we use LMSDiscreteScheduler, let's make sure latents are multiplied by sigmas
if isinstance(self.scheduler, LMSDiscreteScheduler):
latents = latents * self.scheduler.sigmas[0]
Expand All @@ -247,7 +255,7 @@ def __call__(
if accepts_eta:
extra_step_kwargs["eta"] = eta

for i, t in enumerate(self.progress_bar(self.scheduler.timesteps)):
for i, t in enumerate(self.progress_bar(timesteps_tensor)):
# expand the latents if we are doing classifier free guidance
latent_model_input = torch.cat([latents] * 2) if do_classifier_free_guidance else latents
if isinstance(self.scheduler, LMSDiscreteScheduler):
Expand Down Expand Up @@ -278,7 +286,9 @@ def __call__(

# run safety checker
safety_checker_input = self.feature_extractor(self.numpy_to_pil(image), return_tensors="pt").to(self.device)
image, has_nsfw_concept = self.safety_checker(images=image, clip_input=safety_checker_input.pixel_values)
image, has_nsfw_concept = self.safety_checker(
images=image, clip_input=safety_checker_input.pixel_values.to(text_embeddings.dtype)
)

if output_type == "pil":
image = self.numpy_to_pil(image)
Expand Down
Original file line number Diff line number Diff line change
Expand Up @@ -265,7 +265,11 @@ def __call__(
latents = init_latents

t_start = max(num_inference_steps - init_timestep + offset, 0)
for i, t in enumerate(self.progress_bar(self.scheduler.timesteps[t_start:])):
# Some schedulers like PNDM have timesteps as arrays
# It's more optimzed to move all timesteps to correct device beforehand
timesteps_tensor = torch.tensor(self.scheduler.timesteps[t_start:], device=self.device)

for i, t in enumerate(self.progress_bar(timesteps_tensor)):
t_index = t_start + i

# expand the latents if we are doing classifier free guidance
Expand Down
Original file line number Diff line number Diff line change
Expand Up @@ -298,7 +298,11 @@ def __call__(

latents = init_latents
t_start = max(num_inference_steps - init_timestep + offset, 0)
for i, t in tqdm(enumerate(self.scheduler.timesteps[t_start:])):
# Some schedulers like PNDM have timesteps as arrays
# It's more optimzed to move all timesteps to correct device beforehand
timesteps_tensor = torch.tensor(self.scheduler.timesteps[t_start:], device=self.device)

for i, t in tqdm(enumerate(timesteps_tensor)):
t_index = t_start + i
# expand the latents if we are doing classifier free guidance
latent_model_input = torch.cat([latents] * 2) if do_classifier_free_guidance else latents
Expand Down
Loading