-
Notifications
You must be signed in to change notification settings - Fork 169
New issue
Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.
By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.
Already on GitHub? Sign in to your account
Change NF4Tensor dtype and add support for linear #62
Conversation
@@ -546,7 +580,7 @@ class LinearNF4(torch.autograd.Function): | |||
def forward(ctx, input: torch.Tensor, weight: NF4Tensor): | |||
"""Save the quantized nf4 weight for backward pass""" | |||
ctx.nf4_weight = weight | |||
return F.linear(input, weight.get_original_weight()) |
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
@drisspg - So we used to dequantize for each linear call? I guess that makes sense since it's essentially weight only quant.
torchao/dtypes/nf4tensor.py
Outdated
@@ -160,7 +192,7 @@ def __new__( | |||
tensor_meta.original_shape, | |||
tensor_meta.original_strides, | |||
tensor_meta.storage_offset, | |||
dtype=tensor_meta.dtype, | |||
dtype=torch.bits2x4, |
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
for provenance I still dont like this 🙃
I think that nf4tensor's outer wrapper subclass should have the same dtype as the type that it was created from.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
I agree. We need a better extensibility story for dtypes.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
yeah I think we want to deprecate these, why not use torch.uint2?
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
nf4 is a 4bit type. I suppose another mitigation is a type guard at torch_dispatch level and using torch.bits8 just so the allocator will always spit out bytes (not like it has a choice).
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
torch.bits2x4 means 8 bit though, these dtypes (including bits1x8, bits4x2) should be removed actually, since torch.bits8 means the same thing because the meaning is uninterpreted dtypes
so what are you trying to express here? 2 bits * 2 that packed into a 4 bit?
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
this sounds like a uint4 tensor with a different packing format, can you reuse uint4 Tensor as the underlying dtype (by inheriting from UInt4Tensor probably)? can you write down all the use cases for nf4 dtype as well so we get some idea of how we can support it?
bits8 is generally not recommended right now either btw, since all these bit shifting ops etc. are already available in uint8 so we'd recommend uint8 if you want a 8 bit dtype.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
for provenance I still dont like this 🙃. I think that nf4tensor's outer wrapper subclass should have the same dtype as the type that it was created from.
I agree. Having this represent the high precision dtype has worked well for Float8Tensor
.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
Yeah Uint4Tensor is the same as NF4Tensor, AFAIK I think ed copied the packing format from NFTensor in nuggets and that was the basis of uint4tensor.
Nf4Tensor was copied over to ao and not inherited for speed of enabling torchtune. But I agree that NF4 should like inherit from uint4
That being said this same outer tensor dtype issue applies the same for the uint4tensor same as it does this
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
So Float8Tensor's dtype is bfloat16?
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
yes. For float8 tensor specifically, this is required, because we need to trick autograd's assert x.dtype == x.grad.dtype
restriction. But it's also conceptually simple to reason about, "this is an emulation of a bfloat16 tensor with scaled float8".
torchao/dtypes/nf4tensor.py
Outdated
@@ -188,6 +221,7 @@ def __init__( | |||
self.scaler_mean = scaler_mean | |||
self.quantized_data = quantized_data | |||
self.nf4 = nf4 | |||
self.transpose = transpose |
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
If we are doing the lazy transpose you need to update the unlatten and flatten methods
https://github.com/drisspg/transformer_nuggets/pull/24/files#diff-6954806b06a395823616c5a0fefc77b5cf274b3f88a5381c38856cbf74d6f2bfR420
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
Hm, maybe I'll put this into the strides instead and rely on is_contiguous
instead.
inpt_tensor = torch.rand(128, dtype=torch.bfloat16) | ||
inpt_tensor_nf4 = to_nf4(inpt_tensor, 32, 2) | ||
assert type(inpt_tensor_nf4) != torch.Tensor | ||
assert type(inpt_tensor_nf4.to(torch.bfloat16)) == torch.Tensor |
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
Also check the dtype of inpt_tensor_nf4.to(torch.bfloat16)
?
test/dtypes/test_nf4.py
Outdated
out3 = torch.compile(torch.nn.functional.linear, mode='max-autotune')(inp, a_nf4) | ||
|
||
# torch.testing.assert_allclose(out1, out2) | ||
# torch.testing.assert_allclose(out1, out3) |
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
Do these tests pass?
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
Not without custom atol/rtol. But this happens in a test higher up. I need to refactor these tests a bit more before landing.
@@ -428,7 +462,7 @@ def quantize_tensor_nearest( | |||
# pyre-fixme[40]: Static method `dequantize` cannot override a non-static method | |||
# defined in `torch._C.TensorBase`. | |||
def dequantize(value: torch.Tensor, nf4: torch.Tensor) -> torch.Tensor: | |||
"""Dequantize a nf4 value to float16 format""" | |||
"""Dequantize a nf4 value to bfloat16 format""" |
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
Is nf4 tensor still restricted to bf16 only for the higher precision, are there any blockers in supporting fp32?
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
We should be able to support arbitrary precision for conversion, but of course the fidelity of nf4 is independen of the dtype that was passed during construction.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
LG overall from my end
Current behavior
New behavior
The previous behavior happened because the
dtype
of NF4Tensor was bfloat16 and so the conversionto(torch.bfloat16)
was seen as idempotent and didn't actually trigger any underlying functions.