-
Notifications
You must be signed in to change notification settings - Fork 61
Add LLMC integration test #1053
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
Merged
Merged
Changes from all commits
Commits
Show all changes
10 commits
Select commit
Hold shift + click to select a range
a8494ac
add test for llmc
yiliu30 f9e07ca
add llmc
yiliu30 cd16b7d
rm ct
yiliu30 2f8fce1
fix
yiliu30 ce60aee
Merge branch 'main' into add-llmc-test
yiliu30 081926f
merge main
yiliu30 62d9160
fix
yiliu30 1252c26
fix device
yiliu30 0c2ec8d
Merge branch 'main' into add-llmc-test
XuehaoSun 19f5da9
fix
yiliu30 File filter
Filter by extension
Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
There are no files selected for viewing
This file contains hidden or bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
This file contains hidden or bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
This file contains hidden or bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
| Original file line number | Diff line number | Diff line change |
|---|---|---|
| @@ -0,0 +1,93 @@ | ||
| import pytest | ||
| import torch | ||
| from compressed_tensors.quantization import QuantizationArgs, QuantizationScheme | ||
| from llmcompressor import oneshot | ||
| from llmcompressor.modifiers.autoround import AutoRoundModifier | ||
| from transformers import AutoModelForCausalLM, AutoTokenizer | ||
|
|
||
| from auto_round.calib_dataset import get_dataset | ||
|
|
||
| recipe_str = """ | ||
| quant_stage: | ||
| quant_modifiers: | ||
| AutoRoundModifier: | ||
| ignore: ["lm_head"] | ||
| iters: 1 | ||
| config_groups: | ||
| group_0: | ||
| targets: | ||
| - "Linear" | ||
| input_activations: null | ||
| output_activations: null | ||
| weights: | ||
| num_bits: 4 | ||
| type: "int" | ||
| symmetric: true | ||
| strategy: group | ||
| group_size: 128 | ||
| """ | ||
|
|
||
| recipe_modifier_full = AutoRoundModifier( | ||
| ignore=["lm_head"], | ||
| iters=1, | ||
| config_groups={ | ||
| "group_0": QuantizationScheme( | ||
| targets=["Linear"], | ||
| weights=QuantizationArgs(num_bits=4, strategy="group", group_size=128), | ||
| ) | ||
| }, | ||
| ) | ||
|
|
||
|
|
||
| @pytest.mark.parametrize( | ||
| "recipe", | ||
| [ | ||
| recipe_str, | ||
| recipe_modifier_full, | ||
| ], | ||
| ) | ||
| def test_oneshot_application(recipe, tmp_path): | ||
| output = tmp_path / "oneshot_output" | ||
| model = "TinyLlama/TinyLlama-1.1B-Chat-v1.0" | ||
| tokenizer = AutoTokenizer.from_pretrained(model) | ||
| dataset = get_dataset( | ||
| tokenizer=tokenizer, | ||
| seqlen=16, | ||
| nsamples=2, | ||
| ) | ||
| device = "cuda:0" if torch.cuda.is_available() else "cpu" | ||
|
|
||
| oneshot( | ||
| model=model, | ||
| dataset=dataset, | ||
| output_dir=output, | ||
| recipe=recipe, | ||
| ) | ||
| model_loaded = AutoModelForCausalLM.from_pretrained(output, device_map=device) | ||
|
|
||
| # Check that the model is quantized | ||
| # for compression_config - decompress() will attach a quantization_config | ||
| # to the model as we decompress right away | ||
| # for quantization_config - we have CompressedLinear which will only | ||
| # decompress on the forward pass and does not call decompress(). Results | ||
| # in a slightly different parameter tree to access the quant config | ||
| quantization_config = model_loaded.config.quantization_config.quantization_config | ||
| assert quantization_config is not None | ||
|
|
||
| # check config is set properly | ||
| assert "lm_head" in quantization_config.ignore | ||
| assert len(quantization_config.config_groups) == 1 | ||
| quant_scheme = quantization_config.config_groups["group_0"] | ||
| assert isinstance(quant_scheme, QuantizationScheme) | ||
|
|
||
| weight_args = quantization_config.config_groups["group_0"].weights | ||
| assert isinstance(weight_args, QuantizationArgs) | ||
| assert weight_args.num_bits == 4 | ||
|
|
||
| # Check a specific layer is quantized | ||
| targeted_linear_layer = model_loaded.model.layers[2].self_attn.q_proj | ||
| assert hasattr(targeted_linear_layer, "quantization_scheme") | ||
yiliu30 marked this conversation as resolved.
Show resolved
Hide resolved
|
||
|
|
||
| # Check lm-head is not quantized | ||
| not_targeted = model_loaded.lm_head | ||
| assert not hasattr(not_targeted, "quantization_scheme") | ||
Add this suggestion to a batch that can be applied as a single commit.
This suggestion is invalid because no changes were made to the code.
Suggestions cannot be applied while the pull request is closed.
Suggestions cannot be applied while viewing a subset of changes.
Only one suggestion per line can be applied in a batch.
Add this suggestion to a batch that can be applied as a single commit.
Applying suggestions on deleted lines is not supported.
You must change the existing code in this line in order to create a valid suggestion.
Outdated suggestions cannot be applied.
This suggestion has been applied or marked resolved.
Suggestions cannot be applied from pending reviews.
Suggestions cannot be applied on multi-line comments.
Suggestions cannot be applied while the pull request is queued to merge.
Suggestion cannot be applied right now. Please check back later.
Uh oh!
There was an error while loading. Please reload this page.